-
Notifications
You must be signed in to change notification settings - Fork 486
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
Enable eager spmd #7341
Merged
Merged
Enable eager spmd #7341
Changes from all commits
Commits
Show all changes
4 commits
Select commit
Hold shift + click to select a range
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
49 changes: 49 additions & 0 deletions
49
examples/eager/train_decoder_only_eager_spmd_data_parallel.py
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 |
---|---|---|
@@ -0,0 +1,49 @@ | ||
import sys | ||
import os | ||
example_folder = os.path.dirname(os.path.dirname(os.path.abspath(sys.argv[0]))) | ||
sys.path.append(example_folder) | ||
from train_decoder_only_base import TrainDecoderOnlyBase | ||
|
||
import numpy as np | ||
|
||
import torch | ||
import torch_xla | ||
import torch_xla.distributed.spmd as xs | ||
import torch_xla.distributed.parallel_loader as pl | ||
import torch_xla.utils.utils as xu | ||
from torch_xla import runtime as xr | ||
|
||
# Enable the SPMD | ||
xr.use_spmd() | ||
|
||
|
||
# More detailed examaple can be found in https://github.com/pytorch/xla/blob/master/test/spmd/test_train_spmd_imagenet.py | ||
# Check out our user guide in https://github.com/pytorch/xla/blob/master/docs/spmd.md | ||
class TrainDecoderSpmdDDP(TrainDecoderOnlyBase): | ||
|
||
def __init__(self): | ||
super().__init__() | ||
# Shard along batch dimension only | ||
num_devices = xr.global_runtime_device_count() | ||
device_ids = np.arange(num_devices) | ||
mesh_shape = (num_devices,) | ||
mesh = xs.Mesh(device_ids, mesh_shape, ('data',)) | ||
# scale the batch size with num_devices since there will be only one | ||
# process that handles all runtime devices. | ||
self.batch_size *= num_devices | ||
|
||
train_loader = xu.SampleGenerator( | ||
data=(torch.zeros(self.batch_size, self.seq_len, dtype=torch.int64), | ||
torch.zeros(self.batch_size, self.seq_len, dtype=torch.int64)), | ||
sample_count=self.train_dataset_len // self.batch_size) | ||
self.train_device_loader = pl.MpDeviceLoader( | ||
train_loader, | ||
self.device, | ||
# Shard the input's batch dimension along the `data` axis, no sharding along other dimensions | ||
input_sharding=xs.ShardingSpec(mesh, ('data', None))) | ||
|
||
|
||
if __name__ == '__main__': | ||
torch_xla.experimental.eager_mode(True) | ||
spmd_ddp = TrainDecoderSpmdDDP() | ||
spmd_ddp.start_training() |
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 |
---|---|---|
@@ -0,0 +1,44 @@ | ||
import unittest | ||
import sys | ||
|
||
import torch | ||
import torch_xla | ||
import torch_xla.debug.metrics as met | ||
import torch_xla.core.xla_model as xm | ||
import torch_xla.runtime as xr | ||
import torch_xla.distributed.spmd as xs | ||
import numpy as np | ||
|
||
|
||
class Eager(unittest.TestCase): | ||
|
||
@classmethod | ||
def setUpClass(cls): | ||
torch_xla.experimental.eager_mode(True) | ||
xr.use_spmd() | ||
cls.n_devices = xr.global_runtime_device_count() | ||
cls.device_ids = np.array(range(cls.n_devices)) | ||
|
||
def _get_mesh(self, mesh_shape, device_ids=None, axis_names=None): | ||
assert type(mesh_shape) is tuple, 'mesh_shape must be Tuple[int]' | ||
if device_ids is None: | ||
device_ids = self.device_ids | ||
assert len(device_ids) == self.n_devices | ||
return xs.Mesh(device_ids, mesh_shape, axis_names) | ||
|
||
def test_eager_spmd_basic(self): | ||
device = torch_xla.device() | ||
mesh = self._get_mesh((self.n_devices,), axis_names=('data',)) | ||
torch.manual_seed(100) | ||
linear = torch.nn.Linear(10, 20) | ||
input = torch.randn(8, 10) | ||
input_xla = input.to(device) | ||
res = linear(input) | ||
linear.to(device) | ||
res_xla = linear(input_xla) | ||
self.assertTrue(torch.allclose(res, res_xla.cpu(), rtol=1e-3)) | ||
|
||
|
||
if __name__ == '__main__': | ||
test = unittest.main() | ||
sys.exit(0 if test.result.wasSuccessful() else 1) |
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
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
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
Oops, something went wrong.
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.
May I ask why?
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.
It complained about the output tensor already has a sharding and we can't propagate to it. This happens in the backward. I didn't spend enough time to debug it but I don't expect user to actually run eager mode with step fn(forward and backward), I only expect them to run it with some data preprocessing on device so I just quickly unblock myself.