Skip to content

add inputs_embeds to Bart/MBart/Unified_Transformer/Unimo/CodeGen #3769

New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Merged
merged 25 commits into from
Dec 16, 2022

Conversation

Yam0214
Copy link
Contributor

@Yam0214 Yam0214 commented Nov 15, 2022

PR types

New features

PR changes

Models

Description

  1. add inputs_embeds to Bart/MBart/Unified_Transformer/Unimo/CodeGen
  2. force use_cache to False if labels is provided for Bart and MBart Model (in training)

@Yam0214 Yam0214 marked this pull request as draft November 15, 2022 11:21
@Yam0214 Yam0214 marked this pull request as ready for review November 16, 2022 09:23
@Yam0214 Yam0214 changed the title add inputs_embeds to Bart/MBart... add inputs_embeds to Bart/MBart/Unified_Transformer/Unimo/CodeGen Nov 16, 2022
@@ -515,11 +540,13 @@ def set_input_embeddings(self, value):

def forward(
self,
input_ids,
input_ids=None,
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

也修改下文档,加一个 optional,后面的地方也一样。

if attention_mask is None:
assert input_ids is not None, "input_ids should be " "specified when generating attention_mask"
if attention_mask is None and input_ids is not None:
# assert input_ids is not None, "input_ids should be " \
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

这里能否也和 codegen 的修改一样,在 input_ids 为 None 的时候贴一个 warning,然后删除这个注释

elif input_ids is not None:
inputs_sample = input_ids
elif input_embeddings is not None:
inputs_sample = input_embeddings[:, :, -1]
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

这里的 inputs_sample 在其表示意义上有点奇怪,后面的 paddle.expand_as() 能否替换成 paddle.expand(),然后这里改成获取 input_shape,也避免后续反复调用 paddle.shape()

).astype("int64")
else:
logger.warning(
"position_ids or pad_token_ids should be provided when input_embeds is specified, otherwise an unexpected result may be returned"
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

otherwise 这句,说明下是怎样的 position id 吧,an unexpected result 有些笼统了

paddle.cast(input_ids == self.pad_token_id, dtype=paddle.get_default_dtype()).unsqueeze([1, 2])
* -1e4
)
logger.warning("provided inputs_embeds without attention_mask")
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

首字母大写,其次说明下将会使用默认值 None 作为 attention mask,表示不进行 mask 操作。后面的也一样。

elif input_ids is not None:
inputs_sample = input_ids
elif input_embeddings is not None:
inputs_sample = input_embeddings[:, :, -1]
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

这里和上面一样,需要修改下。

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

  1. input_ids 的文档加上 optional
  2. assert input_ids is not None 注释删除,并加上logger
  3. paddle.expand_as 部分替换成 paddle.expand
  4. 补充了 logger.warning的内容,开头大写。

@codecov
Copy link

codecov bot commented Dec 16, 2022

Codecov Report

Merging #3769 (d0aded0) into develop (271f3c1) will increase coverage by 0.10%.
The diff coverage is 72.04%.

@@             Coverage Diff             @@
##           develop    #3769      +/-   ##
===========================================
+ Coverage    32.95%   33.06%   +0.10%     
===========================================
  Files          400      400              
  Lines        56031    56131     +100     
===========================================
+ Hits         18466    18560      +94     
- Misses       37565    37571       +6     
Impacted Files Coverage Δ
paddlenlp/transformers/t5/modeling.py 16.42% <0.00%> (-0.03%) ⬇️
...lenlp/transformers/unified_transformer/modeling.py 81.35% <57.69%> (-1.98%) ⬇️
paddlenlp/transformers/unimo/modeling.py 82.08% <57.69%> (-2.10%) ⬇️
paddlenlp/transformers/mbart/modeling.py 80.73% <73.21%> (-1.26%) ⬇️
paddlenlp/transformers/codegen/modeling.py 88.85% <78.94%> (+<0.01%) ⬆️
paddlenlp/transformers/bart/modeling.py 84.81% <84.21%> (-0.48%) ⬇️
paddlenlp/utils/downloader.py 52.86% <0.00%> (+9.42%) ⬆️

Help us with your feedback. Take ten seconds to tell us how you rate us. Have a feature suggestion? Share it here.

@FrostML FrostML merged commit 1858416 into PaddlePaddle:develop Dec 16, 2022
Sign up for free to join this conversation on GitHub. Already have an account? Sign in to comment
Labels
None yet
Projects
None yet
Development

Successfully merging this pull request may close these issues.

3 participants