-
Notifications
You must be signed in to change notification settings - Fork 2.9k
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
[Dygraph] Support sharding stage2/3+dp in GPT-3 model #2471
Merged
gongel
merged 6 commits into
PaddlePaddle:develop
from
haohongxiang:support_sharding_dp
Sep 29, 2022
Merged
Changes from all commits
Commits
Show all changes
6 commits
Select commit
Hold shift + click to select a range
d13db15
add sharding+dp
haohongxiang 5793563
update
haohongxiang 0025c3d
Merge branch 'develop' into support_sharding_dp
haohongxiang 92910b8
code style check
haohongxiang 2bd8c70
Merge branch 'develop' into support_sharding_dp
1004a03
Merge branch 'develop' into support_sharding_dp
File filter
Filter by extension
Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
There are no files selected for viewing
This file contains bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
Original file line number | Diff line number | Diff line change |
---|---|---|
|
@@ -37,6 +37,8 @@ | |
from paddle.distributed import fleet | ||
from paddle.distributed.fleet.meta_parallel import get_rng_state_tracker | ||
from paddle.distributed.fleet.meta_optimizers.dygraph_optimizer import DygraphShardingOptimizer | ||
from paddle.fluid.dygraph.parallel import sync_params_buffers | ||
from paddle.distributed.fleet.utils.hybrid_parallel_util import fused_allreduce_gradients | ||
|
||
# add sharding stage2/3 | ||
from paddle.distributed.sharding import group_sharded_parallel | ||
|
@@ -151,9 +153,10 @@ def do_train(args): | |
dp_rank = hcg.get_data_parallel_rank() | ||
sharding_rank = hcg.get_sharding_parallel_rank() | ||
|
||
# sharding stage2/3 not support hybrid parallel | ||
# sharding stage2/3 not support hybrid parallel now | ||
if args.sharding_stage in [2, 3]: | ||
assert args.dp_degree == args.mp_degree == args.pp_degree == 1, "sharding stage2/3 will support hybrid parallel later" | ||
assert args.mp_degree == args.pp_degree == 1, "sharding stage2/3 will support tensor/pipeline parallel later" | ||
dp_group = hcg.get_data_parallel_group() | ||
|
||
sharding_size = hcg.get_sharding_parallel_world_size() | ||
data_world_rank = dp_rank * sharding_size + sharding_rank | ||
|
@@ -275,6 +278,11 @@ def do_train(args): | |
# wrap sharding stage2/3 and add collective group | ||
# TODO(Baibaifan): combine ShardingStage1/2/3 and fleet.distributed_model in feature | ||
if args.sharding_stage in [2, 3]: | ||
if args.dp_degree > 1: | ||
sync_params_buffers(model, | ||
comm_group=dp_group, | ||
src_rank=dp_group.ranks[0]) | ||
|
||
scaler = scaler if args.use_pure_fp16 else None | ||
model, optimizer, scaler = wrap_sharding_2_3(model, optimizer, scaler, | ||
args.sharding_offload) | ||
|
@@ -359,6 +367,16 @@ def do_train(args): | |
loss_mbs.backward() | ||
loss = loss + loss_mbs | ||
|
||
if args.sharding_stage in [2, 3] and args.dp_degree > 1: | ||
fused_allreduce_gradients(model.parameters(), hcg) | ||
if args.sharding_stage == 3: | ||
for p in model.parameters(): | ||
if hasattr(p, "bw_storage"): | ||
There was a problem hiding this comment. Choose a reason for hiding this commentThe reason will be displayed to describe this comment to others. Learn more. 这个 |
||
assert p.grad is None, "This case shouldn't happen." | ||
p.bw_storage.scale_(1.0 / dp_group.nranks) | ||
paddle.distributed.all_reduce( | ||
p.bw_storage, group=dp_group) | ||
|
||
if args.use_pure_fp16: | ||
if args.sharding_stage in [2, 3]: | ||
scaler.step(optimizer) | ||
|
Add this suggestion to a batch that can be applied as a single commit.
This suggestion is invalid because no changes were made to the code.
Suggestions cannot be applied while the pull request is closed.
Suggestions cannot be applied while viewing a subset of changes.
Only one suggestion per line can be applied in a batch.
Add this suggestion to a batch that can be applied as a single commit.
Applying suggestions on deleted lines is not supported.
You must change the existing code in this line in order to create a valid suggestion.
Outdated suggestions cannot be applied.
This suggestion has been applied or marked resolved.
Suggestions cannot be applied from pending reviews.
Suggestions cannot be applied on multi-line comments.
Suggestions cannot be applied while the pull request is queued to merge.
Suggestion cannot be applied right now. Please check back later.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
这里
fused_allreduce_gradient
和下面都加一下注释吧。不清楚是做了啥通信