Skip to content
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

support sharding in fp16 on xpu, #48897

Merged
merged 3 commits into from
Dec 12, 2022
Merged
Show file tree
Hide file tree
Changes from all commits
Commits
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
Original file line number Diff line number Diff line change
Expand Up @@ -220,7 +220,8 @@ def unscale_method(self, optimizer):
temp_found_inf_fp16 = to_variable(np.array([0]).astype(np.bool_))
temp_found_inf_fp32 = to_variable(np.array([0]).astype(np.bool_))

device = "cpu" if optimizer.offload else "gpu"
device = paddle.get_device().split(":")[0]
device = "cpu" if optimizer.offload else device
dev_id = (
0 if device == "cpu" else int(paddle.get_device().split(":")[1])
)
Expand All @@ -245,8 +246,9 @@ def unscale_method(self, optimizer):
is_found_inf = paddle.to_tensor([self._found_inf], dtype="int32")

paddle.distributed.all_reduce(
is_found_inf, op=paddle.distributed.ReduceOp.MAX, group=None
is_found_inf, op=paddle.distributed.ReduceOp.SUM, group=None
)

self._found_inf = is_found_inf.numpy()[0]

scaler._unscale = MethodType(unscale_method, scaler)
Expand Down
Original file line number Diff line number Diff line change
Expand Up @@ -344,7 +344,7 @@ def nan_inf(self):
scaled_loss = scaler.scale(loss)
scaled_loss.backward()
optimize_ops, params_grads = scaler.minimize(optimizer, scaled_loss)
self.assertEqual(scaler._found_inf.numpy() == 1, True)
self.assertEqual(scaler._found_inf.numpy() >= 1, True)

for param in model.parameters():
# param not update when tensor contains nan or inf
Expand Down
Original file line number Diff line number Diff line change
Expand Up @@ -343,7 +343,7 @@ def nan_inf(self):
scaled_loss = scaler.scale(loss)
scaled_loss.backward()
optimize_ops, params_grads = scaler.minimize(optimizer, scaled_loss)
self.assertEqual(scaler._found_inf.numpy() == 1, True)
self.assertEqual(scaler._found_inf.numpy() >= 1, True)

for param in model.parameters():
# param not update when tensor contains nan or inf
Expand Down