-
Notifications
You must be signed in to change notification settings - Fork 3.3k
/
test_fsdp.py
587 lines (462 loc) · 23 KB
/
test_fsdp.py
1
2
3
4
5
6
7
8
9
10
11
12
13
14
15
16
17
18
19
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
55
56
57
58
59
60
61
62
63
64
65
66
67
68
69
70
71
72
73
74
75
76
77
78
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
104
105
106
107
108
109
110
111
112
113
114
115
116
117
118
119
120
121
122
123
124
125
126
127
128
129
130
131
132
133
134
135
136
137
138
139
140
141
142
143
144
145
146
147
148
149
150
151
152
153
154
155
156
157
158
159
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
177
178
179
180
181
182
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
220
221
222
223
224
225
226
227
228
229
230
231
232
233
234
235
236
237
238
239
240
241
242
243
244
245
246
247
248
249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
265
266
267
268
269
270
271
272
273
274
275
276
277
278
279
280
281
282
283
284
285
286
287
288
289
290
291
292
293
294
295
296
297
298
299
300
301
302
303
304
305
306
307
308
309
310
311
312
313
314
315
316
317
318
319
320
321
322
323
324
325
326
327
328
329
330
331
332
333
334
335
336
337
338
339
340
341
342
343
344
345
346
347
348
349
350
351
352
353
354
355
356
357
358
359
360
361
362
363
364
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
385
386
387
388
389
390
391
392
393
394
395
396
397
398
399
400
401
402
403
404
405
406
407
408
409
410
411
412
413
414
415
416
417
418
419
420
421
422
423
424
425
426
427
428
429
430
431
432
433
434
435
436
437
438
439
440
441
442
443
444
445
446
447
448
449
450
451
452
453
454
455
456
457
458
459
460
461
462
463
464
465
466
467
468
469
470
471
472
473
474
475
476
477
478
479
480
481
482
483
484
485
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
501
502
503
504
505
506
507
508
509
510
511
512
513
514
515
516
517
518
519
520
521
522
523
524
525
526
527
528
529
530
531
532
533
534
535
536
537
538
539
540
541
542
543
544
545
546
547
548
549
550
551
552
553
554
555
556
557
558
559
560
561
562
563
564
565
566
567
568
569
570
571
572
573
574
575
576
577
578
579
580
581
582
583
584
585
586
587
import os
from contextlib import nullcontext
from datetime import timedelta
from functools import partial
from typing import Any, Callable, Dict, Optional
from unittest import mock
from unittest.mock import ANY, Mock
import pytest
import torch
import torch.nn as nn
from lightning.fabric.plugins.environments import LightningEnvironment
from lightning.fabric.utilities.imports import _TORCH_GREATER_EQUAL_1_12, _TORCH_GREATER_EQUAL_2_0
from lightning.pytorch import Trainer
from lightning.pytorch.callbacks import ModelCheckpoint
from lightning.pytorch.demos.boring_classes import BoringModel
from lightning.pytorch.plugins.precision.fsdp import FSDPMixedPrecisionPlugin
from lightning.pytorch.strategies import FSDPStrategy
from lightning.pytorch.utilities.exceptions import MisconfigurationException
from tests_pytorch.helpers.runif import RunIf
if _TORCH_GREATER_EQUAL_1_12:
from torch.distributed.fsdp.fully_sharded_data_parallel import CPUOffload, FullyShardedDataParallel, MixedPrecision
from torch.distributed.fsdp.wrap import size_based_auto_wrap_policy, wrap
else:
size_based_auto_wrap_policy = object
if _TORCH_GREATER_EQUAL_2_0:
from torch.distributed.fsdp.wrap import _FSDPPolicy
else:
_FSDPPolicy = object
class TestFSDPModel(BoringModel):
def __init__(self):
super().__init__()
self.layer: Optional[torch.nn.Module] = None
def _init_model(self) -> None:
self.layer = torch.nn.Sequential(torch.nn.Linear(32, 32), torch.nn.ReLU(), torch.nn.Linear(32, 2))
def setup(self, stage: str) -> None:
if self.layer is None:
self._init_model()
def configure_sharded_model(self) -> None:
# the model is already wrapped with FSDP: no need to wrap again!
if isinstance(self.layer, FullyShardedDataParallel):
return
for i, layer in enumerate(self.layer):
if i % 2 == 0:
self.layer[i] = wrap(layer)
self.layer = wrap(self.layer)
def on_load_checkpoint(self, checkpoint: Dict[str, Any]) -> None:
# when loading full state dict, we first need to create a new unwrapped model
self._init_model()
def configure_optimizers(self):
# There is some issue with SGD optimizer state in FSDP
return torch.optim.AdamW(self.layer.parameters(), lr=0.1)
def on_train_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def on_test_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def on_validation_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def on_predict_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def _assert_layer_fsdp_instance(self) -> None:
assert isinstance(self.layer, FullyShardedDataParallel)
assert isinstance(self.trainer.strategy.precision_plugin, FSDPMixedPrecisionPlugin)
if self.trainer.precision == "16-mixed":
param_dtype = torch.float32
reduce_dtype = buffer_dtype = torch.float16
elif self.trainer.precision == "bf16-mixed":
param_dtype = torch.float32
reduce_dtype = buffer_dtype = torch.bfloat16
elif self.trainer.precision == "16-true":
param_dtype = reduce_dtype = buffer_dtype = torch.float16
elif self.trainer.precision == "bf16-true":
param_dtype = reduce_dtype = buffer_dtype = torch.bfloat16
else:
raise ValueError(f"Unknown precision {self.trainer.precision}")
assert self.layer.mixed_precision.param_dtype == param_dtype
assert self.layer.mixed_precision.reduce_dtype == reduce_dtype
assert self.layer.mixed_precision.buffer_dtype == buffer_dtype
for layer_num in [0, 2]:
assert isinstance(self.layer.module[layer_num], FullyShardedDataParallel)
assert self.layer[layer_num].mixed_precision.param_dtype == param_dtype
assert self.layer[layer_num].mixed_precision.reduce_dtype == reduce_dtype
assert self.layer[layer_num].mixed_precision.buffer_dtype == buffer_dtype
class TestBoringModel(BoringModel):
def __init__(self, wrap_min_params: int = 2, automatic_optimization: bool = True):
super().__init__()
self.save_hyperparameters()
self.layer = torch.nn.Sequential(torch.nn.Linear(32, 32), torch.nn.ReLU(), torch.nn.Linear(32, 2))
self.should_be_wrapped = [(32 * 32 + 32) > wrap_min_params, None, (32 * 2 + 2) > wrap_min_params]
self.automatic_optimization = automatic_optimization
def configure_optimizers(self):
parameters = self.parameters() if _TORCH_GREATER_EQUAL_2_0 else self.trainer.model.parameters()
# There are some issue when we are trying to store SGD optimizer state in FSDP
return torch.optim.AdamW(parameters, lr=0.1)
class TestFSDPModelAutoWrapped(TestBoringModel):
def on_train_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def on_test_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def on_validation_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def on_predict_batch_end(self, *_) -> None:
self._assert_layer_fsdp_instance()
def _assert_layer_fsdp_instance(self) -> None:
assert isinstance(self.layer, torch.nn.Sequential)
assert isinstance(self.trainer.strategy.precision_plugin, FSDPMixedPrecisionPlugin)
if self.trainer.precision == "16-mixed":
param_dtype = torch.float32
reduce_dtype = buffer_dtype = torch.float16
elif self.trainer.precision == "bf16-mixed":
param_dtype = torch.float32
reduce_dtype = buffer_dtype = torch.bfloat16
elif self.trainer.precision == "16-true":
param_dtype = reduce_dtype = buffer_dtype = torch.float16
elif self.trainer.precision == "bf16-true":
param_dtype = reduce_dtype = buffer_dtype = torch.bfloat16
else:
raise ValueError(f"Unknown precision {self.trainer.precision}")
for layer_num in [0, 2]:
if not self.should_be_wrapped[layer_num]:
# this layer is not wrapped
assert not isinstance(self.layer[layer_num], FullyShardedDataParallel)
continue
assert isinstance(self.layer[layer_num], FullyShardedDataParallel)
assert self.layer[layer_num].mixed_precision.param_dtype == param_dtype
assert self.layer[layer_num].mixed_precision.reduce_dtype == reduce_dtype
assert self.layer[layer_num].mixed_precision.buffer_dtype == buffer_dtype
def _run_multiple_stages(trainer, model, model_path: Optional[str] = None):
trainer.fit(model)
model_path = trainer.strategy.broadcast(model_path)
model_path = model_path if model_path else trainer.checkpoint_callback.last_model_path
trainer.save_checkpoint(model_path, weights_only=True)
_assert_save_equality(trainer, model_path, cls=model.__class__)
with torch.inference_mode():
# Test entry point
trainer.test(model) # model is wrapped, will not call `configure_sharded_model`
# provide model path, will create a new unwrapped model and load and then call `configure_shared_model` to wrap
trainer.test(ckpt_path=model_path)
# Predict entry point
trainer.predict(model) # model is wrapped, will not call `configure_sharded_model`
# provide model path, will create a new unwrapped model and load and then call `configure_shared_model` to wrap
trainer.predict(ckpt_path=model_path)
def _assert_save_equality(trainer, ckpt_path, cls=TestFSDPModel):
# Use FullySharded to get the state dict for the sake of comparison
model_state_dict = trainer.strategy.lightning_module_state_dict()
if trainer.is_global_zero:
saved_model = cls.load_from_checkpoint(ckpt_path)
# Assert model parameters are identical after loading
for ddp_param, shard_param in zip(model_state_dict.values(), saved_model.state_dict().values()):
assert torch.equal(ddp_param, shard_param)
@RunIf(min_torch="1.12")
def test_invalid_on_cpu(tmpdir):
"""Test to ensure that we raise Misconfiguration for FSDP on CPU."""
with pytest.raises(
MisconfigurationException,
match=f"You selected strategy to be `{FSDPStrategy.strategy_name}`, but GPU accelerator is not used.",
):
trainer = Trainer(default_root_dir=tmpdir, fast_dev_run=True, strategy="fsdp")
assert isinstance(trainer.strategy, FSDPStrategy)
trainer.strategy.setup_environment()
@RunIf(min_torch="1.12", min_cuda_gpus=1)
@pytest.mark.parametrize(
("precision", "expected"),
[
("16-mixed", (torch.float32, torch.float16, torch.float16)),
("bf16-mixed", (torch.float32, torch.bfloat16, torch.bfloat16)),
# TODO: add 16-true and bf16-true once supported
],
)
def test_precision_plugin_config(precision, expected):
plugin = FSDPMixedPrecisionPlugin(precision=precision, device="cuda")
config = plugin.mixed_precision_config
assert config.param_dtype == expected[0]
assert config.buffer_dtype == expected[1]
assert config.reduce_dtype == expected[2]
@RunIf(min_torch="1.12")
def test_fsdp_custom_mixed_precision(tmpdir):
"""Test to ensure that passing a custom mixed precision config works."""
config = MixedPrecision()
strategy = FSDPStrategy(mixed_precision=config)
assert strategy.mixed_precision_config == config
@RunIf(min_cuda_gpus=2, skip_windows=True, standalone=True, min_torch="1.12")
def test_fsdp_strategy_sync_batchnorm(tmpdir):
"""Test to ensure that sync_batchnorm works when using FSDP and GPU, and all stages can be run."""
model = TestFSDPModel()
trainer = Trainer(
default_root_dir=tmpdir,
accelerator="gpu",
devices=2,
strategy="fsdp",
precision="16-mixed",
max_epochs=1,
sync_batchnorm=True,
)
_run_multiple_stages(trainer, model, os.path.join(tmpdir, "last.ckpt"))
@RunIf(min_cuda_gpus=1, skip_windows=True, standalone=True, min_torch="1.12")
@pytest.mark.parametrize("precision", ["16-mixed", pytest.param("bf16-mixed", marks=RunIf(bf16_cuda=True))])
def test_fsdp_strategy_checkpoint(tmpdir, precision):
"""Test to ensure that checkpoint is saved correctly when using a single GPU, and all stages can be run."""
model = TestFSDPModel()
trainer = Trainer(
default_root_dir=tmpdir, accelerator="gpu", devices=1, strategy="fsdp", precision=precision, max_epochs=1
)
_run_multiple_stages(trainer, model, os.path.join(tmpdir, "last.ckpt"))
class CustomWrapPolicy(_FSDPPolicy):
"""This is a wrapper around :func:`_module_wrap_policy`."""
def __init__(self, min_num_params: int):
self._policy: Callable = partial(size_based_auto_wrap_policy, min_num_params=min_num_params)
@property
def policy(self):
return self._policy
custom_fsdp_policy = CustomWrapPolicy(min_num_params=2)
if _TORCH_GREATER_EQUAL_2_0:
def custom_auto_wrap_policy(
module,
recurse,
nonwrapped_numel: int,
) -> bool:
return nonwrapped_numel >= 2
else:
def custom_auto_wrap_policy(
module,
recurse,
unwrapped_params: int,
) -> bool:
return unwrapped_params >= 2
@RunIf(min_cuda_gpus=2, skip_windows=True, standalone=True, min_torch="1.12")
@pytest.mark.parametrize("wrap_min_params", [2, 1024, 100000000])
def test_fsdp_strategy_full_state_dict(tmpdir, wrap_min_params):
"""Test to ensure that the full state dict is extracted when using FSDP strategy.
Based on `wrap_min_params`, the model will be fully wrapped, half wrapped, and not wrapped at all.
"""
model = TestFSDPModelAutoWrapped(wrap_min_params=wrap_min_params)
correct_state_dict = model.state_dict() # State dict before wrapping
strategy = FSDPStrategy(auto_wrap_policy=partial(size_based_auto_wrap_policy, min_num_params=wrap_min_params))
trainer = Trainer(
default_root_dir=tmpdir, accelerator="gpu", devices=2, strategy=strategy, precision="16-mixed", max_epochs=1
)
trainer.fit(model)
full_state_dict = trainer.strategy.lightning_module_state_dict()
if trainer.global_rank != 0:
assert len(full_state_dict) == 0
return
# State dict should contain same number of keys
assert len(correct_state_dict) == len(full_state_dict)
# OrderedDict should return the same keys in the same order
assert all(_ex == _co for _ex, _co in zip(full_state_dict.keys(), correct_state_dict.keys()))
@RunIf(min_cuda_gpus=2, skip_windows=True, standalone=True, min_torch="1.12")
@pytest.mark.parametrize(
("model", "strategy", "strategy_cfg"),
[
pytest.param(TestFSDPModel(), "fsdp", None, id="manually_wrapped"),
pytest.param(
TestFSDPModelAutoWrapped(),
FSDPStrategy,
{"auto_wrap_policy": custom_auto_wrap_policy},
marks=RunIf(max_torch="2.0.0"),
id="autowrap_1x",
),
pytest.param(
TestFSDPModelAutoWrapped(),
FSDPStrategy,
{"auto_wrap_policy": custom_auto_wrap_policy},
marks=RunIf(min_torch="2.0.0"),
id="autowrap_2x",
),
pytest.param(
TestFSDPModelAutoWrapped(),
FSDPStrategy,
{"auto_wrap_policy": custom_fsdp_policy, "use_orig_params": True},
marks=RunIf(min_torch="2.0.0"),
id="autowrap_use_orig_params",
),
],
)
def test_fsdp_checkpoint_multi_gpus(tmpdir, model, strategy, strategy_cfg):
"""Test to ensure that checkpoint is saved correctly when using multiple GPUs, and all stages can be run."""
ck = ModelCheckpoint(save_last=True)
strategy_cfg = strategy_cfg or {}
if not isinstance(strategy, str):
strategy = strategy(**strategy_cfg)
trainer = Trainer(
default_root_dir=tmpdir,
accelerator="gpu",
devices=2,
strategy=strategy,
precision="16-mixed",
max_epochs=1,
limit_train_batches=2,
limit_val_batches=2,
limit_test_batches=2,
limit_predict_batches=2,
callbacks=[ck],
)
_run_multiple_stages(trainer, model)
@RunIf(min_cuda_gpus=1, skip_windows=True, standalone=True, min_torch="1.12")
def test_invalid_parameters_in_optimizer():
trainer = Trainer(
strategy="fsdp",
accelerator="cuda",
devices=1,
fast_dev_run=1,
)
error_context = (
nullcontext()
if _TORCH_GREATER_EQUAL_2_0
else pytest.raises(ValueError, match="The optimizer does not seem to reference any FSDP parameters")
)
class EmptyParametersModel(BoringModel):
def configure_optimizers(self):
return torch.optim.Adam(self.parameters(), lr=1e-2)
model = EmptyParametersModel()
with error_context:
trainer.fit(model)
class NoFlatParametersModel(BoringModel):
def configure_optimizers(self):
layer = torch.nn.Linear(4, 5)
return torch.optim.Adam(layer.parameters(), lr=1e-2)
model = NoFlatParametersModel()
with error_context:
trainer.fit(model)
@RunIf(min_torch="1.12")
@mock.patch("lightning.pytorch.strategies.fsdp._TORCH_GREATER_EQUAL_1_13", False)
def test_fsdp_activation_checkpointing_support():
"""Test that we error out if activation checkpointing requires a newer PyTorch version."""
with pytest.raises(ValueError, match="Activation checkpointing requires torch >= 1.13.0"):
FSDPStrategy(activation_checkpointing=Mock())
@RunIf(min_torch="1.13")
def test_fsdp_activation_checkpointing():
"""Test that the FSDP strategy can apply activation checkpointing to the given layers."""
class Block1(nn.Linear):
pass
class Block2(nn.Linear):
pass
class Model(BoringModel):
def __init__(self):
super().__init__()
self.layer0 = nn.Sequential(Block1(4, 4), Block1(5, 5))
self.layer1 = Block2(2, 2)
self.layer2 = nn.Linear(3, 3)
strategy = FSDPStrategy(activation_checkpointing=Block1)
assert strategy._activation_checkpointing == [Block1]
strategy = FSDPStrategy(activation_checkpointing=[Block1, Block2])
assert strategy._activation_checkpointing == [Block1, Block2]
model = Model()
strategy._parallel_devices = [torch.device("cuda", 0)]
strategy._lightning_module = model
strategy._process_group = Mock()
with mock.patch("lightning.pytorch.strategies.fsdp.FullyShardedDataParallel") as fsdp_mock, mock.patch(
"torch.distributed.algorithms._checkpoint.checkpoint_wrapper.apply_activation_checkpointing"
) as ckpt_mock:
strategy._setup_model(model)
ckpt_mock.assert_called_with(fsdp_mock(), checkpoint_wrapper_fn=ANY, check_fn=ANY)
@RunIf(min_torch="1.12")
def test_fsdp_strategy_cpu_offload():
"""Test the different ways cpu offloading can be enabled."""
# bool
strategy = FSDPStrategy(cpu_offload=True)
assert strategy.cpu_offload == CPUOffload(offload_params=True)
# dataclass
config = CPUOffload()
strategy = FSDPStrategy(cpu_offload=config)
assert strategy.cpu_offload == config
@RunIf(min_torch="1.12")
def test_fsdp_use_orig_params():
"""Test that Lightning enables `use_orig_params` in PyTorch >= 2.0."""
with mock.patch("lightning.pytorch.strategies.fsdp._TORCH_GREATER_EQUAL_2_0", False):
strategy = FSDPStrategy()
assert "use_orig_params" not in strategy.kwargs
with mock.patch("lightning.pytorch.strategies.fsdp._TORCH_GREATER_EQUAL_2_0", True):
strategy = FSDPStrategy()
assert strategy.kwargs["use_orig_params"]
strategy = FSDPStrategy(use_orig_params=False)
assert not strategy.kwargs["use_orig_params"]
@RunIf(min_torch="1.12")
@mock.patch("torch.distributed.init_process_group")
def test_set_timeout(init_process_group_mock):
"""Test that the timeout gets passed to the ``torch.distributed.init_process_group`` function."""
test_timedelta = timedelta(seconds=30)
strategy = FSDPStrategy(timeout=test_timedelta, parallel_devices=[torch.device("cpu")])
strategy.cluster_environment = LightningEnvironment()
strategy.accelerator = Mock()
strategy.setup_environment()
process_group_backend = strategy._get_process_group_backend()
global_rank = strategy.cluster_environment.global_rank()
world_size = strategy.cluster_environment.world_size()
init_process_group_mock.assert_called_with(
process_group_backend, rank=global_rank, world_size=world_size, timeout=test_timedelta
)
@RunIf(min_cuda_gpus=2, skip_windows=True, standalone=True, min_torch="1.12")
@pytest.mark.parametrize("wrap_min_params", [2, 1024, 100000000])
def test_fsdp_strategy_save_optimizer_states(tmpdir, wrap_min_params):
"""Test to ensure that the full state dict and optimizer states is saved when using FSDP strategy.
Based on `wrap_min_params`, the model will be fully wrapped, half wrapped, and not wrapped at all. If the model can
be restored to DDP, it means that the optimizer states were saved correctly.
"""
model = TestFSDPModelAutoWrapped(wrap_min_params=wrap_min_params)
strategy = FSDPStrategy(auto_wrap_policy=partial(size_based_auto_wrap_policy, min_num_params=wrap_min_params))
trainer = Trainer(
default_root_dir=tmpdir, accelerator="gpu", devices=2, strategy=strategy, precision="16-mixed", max_epochs=1
)
trainer.fit(model)
model_path = os.path.join(tmpdir, "last.ckpt")
model_path = trainer.strategy.broadcast(model_path)
trainer.save_checkpoint(model_path)
model_state_dict = trainer.strategy.lightning_module_state_dict()
optimizer_state_dict = trainer.strategy.optimizer_state(model.optimizers())
if trainer.global_rank != 0:
assert len(model_state_dict) == 0
# restore model to ddp, disable automatic_optimization to avoid optimizer state / model state mismatch
model = TestBoringModel(automatic_optimization=False)
trainer = Trainer(
default_root_dir=tmpdir, accelerator="gpu", devices=2, strategy="ddp", precision="16-mixed", max_epochs=1
)
# This step will restore the model and optimizer states
trainer.fit(model, ckpt_path=model_path)
# Get the model and optimizer states from the restored ddp model
restored_model_state_dict = trainer.strategy.lightning_module_state_dict()
restored_optimizer_state_dict = trainer.strategy.optimizer_state(model.optimizers())
if trainer.global_rank != 0:
return
# assert everything is the same
assert len(model_state_dict) == len(restored_model_state_dict)
assert len(optimizer_state_dict) == len(restored_optimizer_state_dict)
torch.testing.assert_close(model_state_dict, restored_model_state_dict, atol=0, rtol=0)
torch.testing.assert_close(optimizer_state_dict, restored_optimizer_state_dict, atol=0, rtol=0)
@RunIf(min_cuda_gpus=2, skip_windows=True, standalone=True, min_torch="1.12")
@pytest.mark.parametrize("wrap_min_params", [2, 1024, 100000000])
def test_fsdp_strategy_load_optimizer_states(tmpdir, wrap_min_params):
"""Test to ensure that the full state dict and optimizer states can be load when using FSDP strategy.
Based on `wrap_min_params`, the model will be fully wrapped, half wrapped, and not wrapped at all. If the DDP model
can be restored to FSDP, it means that the optimizer states were restored correctly.
"""
# restore model to ddp, disable automatic_optimization to avoid optimizer state / model state mismatch
model = TestBoringModel()
trainer = Trainer(
default_root_dir=tmpdir, accelerator="gpu", devices=2, strategy="ddp", precision="16-mixed", max_epochs=1
)
# This step will restore the model and optimizer states
trainer.fit(model)
model_path = os.path.join(tmpdir, "last.ckpt")
model_path = trainer.strategy.broadcast(model_path)
trainer.save_checkpoint(model_path)
# Get the model and optimizer states from the restored ddp model
model_state_dict = trainer.strategy.lightning_module_state_dict()
optimizer_state_dict = trainer.strategy.optimizer_state(model.optimizers())
# Build a new FSDP model, without automatic_optimization
model = TestFSDPModelAutoWrapped(wrap_min_params=wrap_min_params, automatic_optimization=False)
strategy = FSDPStrategy(auto_wrap_policy=partial(size_based_auto_wrap_policy, min_num_params=wrap_min_params))
trainer = Trainer(
default_root_dir=tmpdir, accelerator="gpu", devices=2, strategy=strategy, precision="16-mixed", max_epochs=1
)
trainer.fit(model, ckpt_path=model_path)
restored_model_state_dict = trainer.strategy.lightning_module_state_dict()
restored_optimizer_state_dict = trainer.strategy.optimizer_state(model.optimizers())
if trainer.global_rank != 0:
assert len(restored_model_state_dict) == 0
return
# assert everything is the same
assert len(model_state_dict) == len(restored_model_state_dict)
assert len(optimizer_state_dict) == len(restored_optimizer_state_dict)
torch.testing.assert_close(model_state_dict, restored_model_state_dict, atol=0, rtol=0)
torch.testing.assert_close(optimizer_state_dict, restored_optimizer_state_dict, atol=0, rtol=0)