Skip to content

Commit b5e09d0

Browse files
committedFeb 12, 2025
Integrate TorchFT
**Summary** This is a WIP TorchFT integration PR. **Current Issues** This doesn't work at this moment as there are hanged groups when a new group joins. **Issue 1:** ~Group 0 and group 1 will hang during the first `should_commit` after group 1 applying the pending state_dict from group 0.~ Fixed with: pytorch/torchft#83 **Issue 2:** ~Group 0 and group 1 will pass the `should_commit` but group 0 needs healing which is wrong and the healing process will cause another hang.~ Fixed with: pytorch/torchft#83 **Issue 3:** ~The byproduct of issue 1 and issue 2: group 1 will continue to print out~ ``` [rank0]:devgpu051:76838:80357 [0] misc/socket.cc:50 NCCL WARN socketProgress: Connection closed by remote peer devgpu051.cln3.svc.fbinfra.net<33618> ``` Fixed with pytorch/torchft#91 and several other fixes. **Issue 4:** When there are 3 groups, everyone requests the state dict every step. ***How to reproduce?*** Using the `Reproduce steps` to run 2 groups, then add another group by modifying the command. Seems to be fixed, will need more tests. **Issue 5:** Hang will happen if using functional collective. ***How to reproduce?*** Pull the latest version of this PR and comment out line 41 and uncomment line 42 in `torchtitan/utils.py` **Reproduce steps:** 1. Patch TorchFT with pytorch/torchft#82 2. Execute lighthouse 3. Execute the following command in one terminal: ``` TORCHFT_MANAGER_PORT=29520 REPLICA_GROUP_ID=0 CUDA_VISIBLE_DEVICES=0,1 NGPU=2 ./run_llama_train.sh --training.data_parallel_shard_degree=2 --experimental.enable_torchft --experimental.ft_replica_group_id=0 ``` 4. Wait 10 seconds, execute following command in another terminal: ``` TORCHFT_MANAGER_PORT=29522 REPLICA_GROUP_ID=1 CUDA_VISIBLE_DEVICES=2,3 NGPU=2 ./run_llama_train.sh --training.data_parallel_shard_degree=2 --experimental.enable_torchft --experimental.ft_replica_group_id=1 ``` ghstack-source-id: 088581cceee2c523f2a4ea358f334a0b1cce3927 Pull Request resolved: #834
1 parent 58ab4d1 commit b5e09d0

File tree

8 files changed

+269
-63
lines changed

8 files changed

+269
-63
lines changed
 

‎run_llama_train.sh

+4
Original file line numberDiff line numberDiff line change
@@ -19,7 +19,11 @@ if [ $# -ne 0 ]; then
1919
overrides="$*"
2020
fi
2121

22+
TORCHFT_MANAGER_PORT=${TORCHFT_MANAGER_PORT:-"29512"}
23+
2224
PYTORCH_CUDA_ALLOC_CONF="expandable_segments:True" \
25+
TORCHFT_LIGHTHOUSE=http://localhost:29510 \
26+
TORCHFT_MANAGER_PORT=${TORCHFT_MANAGER_PORT} \
2327
torchrun --nproc_per_node=${NGPU} --rdzv_backend c10d --rdzv_endpoint="localhost:0" \
2428
--local-ranks-filter ${LOG_RANK} --role rank --tee 3 \
2529
train.py --job.config_file ${CONFIG_FILE} $overrides

‎torchtitan/checkpoint.py

+117-50
Original file line numberDiff line numberDiff line change
@@ -19,6 +19,7 @@
1919
import torch.distributed as dist
2020
import torch.distributed.checkpoint as dcp
2121
import torch.nn as nn
22+
from torch.distributed._state_dict_utils import _copy_state_dict, _create_cpu_state_dict
2223
from torch.distributed.checkpoint.state_dict import (
2324
get_model_state_dict,
2425
set_model_state_dict,
@@ -144,13 +145,18 @@ def __init__(
144145
lr_schedulers: LRSchedulersContainer,
145146
states: Dict[str, Any],
146147
job_config: JobConfig,
148+
ft_manager: Optional[Any] = None,
147149
) -> None:
148150
ckpt_config = job_config.checkpoint
149151
self.enable_checkpoint = ckpt_config.enable_checkpoint
150-
self.keep_latest_k = ckpt_config.keep_latest_k
152+
self.ft_manager = ft_manager
153+
self.enable_staging = (
154+
self.enable_checkpoint and async_mode == AsyncMode.ASYNC_WITH_PINNED_MEM
155+
) or self.ft_manager
151156

152-
if not self.enable_checkpoint:
157+
if not self.enable_checkpoint and self.ft_manager is None:
153158
return
159+
154160
"""
155161
Note: Pipeline Parallelism and Virtual Stages
156162
@@ -185,6 +191,13 @@ def __init__(
185191
}
186192
)
187193

194+
async_mode = ckpt_config.async_mode.lower()
195+
self.staging = False
196+
self.sending_to_checkpoint_mp = False
197+
self.staging_id = None
198+
self.cpu_offload_state_dict = None
199+
self.staging_stream = torch.cuda.Stream() if self.enable_staging else None
200+
188201
self.folder = os.path.join(job_config.job.dump_folder, ckpt_config.folder)
189202
self.interval_type = (
190203
IntervalType.SECONDS
@@ -199,6 +212,7 @@ def __init__(
199212
if async_mode == AsyncMode.ASYNC or self.interval_type == IntervalType.SECONDS:
200213
self.pg = dist.new_group(backend="gloo")
201214

215+
self.keep_latest_k = ckpt_config.keep_latest_k
202216
self.model_weights_only = ckpt_config.model_weights_only
203217
self.export_dtype = TORCH_DTYPE_MAP[ckpt_config.export_dtype]
204218
self.exclude_from_loading = ckpt_config.exclude_from_loading
@@ -223,10 +237,6 @@ def __init__(
223237
daemon=True,
224238
)
225239
self.mp.start()
226-
self.cpu_offload_state_dict = None
227-
self.staging = False
228-
self.staging_id = None
229-
self.staging_stream = torch.cuda.Stream()
230240
else:
231241
raise ValueError(f"Unkown checkpoint async_mode {ckpt_config.async_mode}")
232242

@@ -240,8 +250,61 @@ def __del__(self):
240250
self.mp.join()
241251

242252
def reset(self) -> None:
253+
# We need to stage the local state if another replicate joins during the
254+
# first step.
255+
if self.ft_manager:
256+
self.cpu_staging(None)
243257
self.begin_time = time.monotonic()
244258

259+
def _initialize_states(
260+
self,
261+
states: Dict[str, Any],
262+
dataloader: DataLoader,
263+
model_parts: List[nn.Module],
264+
optimizers: OptimizersContainer,
265+
lr_schedulers: LRSchedulersContainer,
266+
) -> None:
267+
"""
268+
Note: Pipeline Parallelism and Virtual Stages
269+
270+
1. Even for simple PP schedules, there is a separate optimizer each PP rank.
271+
rank0's optimizer would have a param_group[0] which refers to layers.0 in the
272+
original model. rank1's would _also_ have a param_group[0], since it's index based,
273+
but referring to layers.1.
274+
When saving, these collide and one of them is lost. Then when reloading, only one
275+
stage can restore its optimizer states, others will error.
276+
277+
The solution to this problem is optimizer flattening: it landed in #127071
278+
and is enabled in TorchTitan by passing the 'flatten_optimizer_state_dict'
279+
kwarg to DCP functions called in the OptimizerContainer.
280+
281+
2. With complex PP schedules, we have multiple model chunks per pp rank. This
282+
compounds challenge (1) by also requiring us to reason about multiple 'optim'
283+
objects locally.
284+
285+
We solve this in the Model and Optimizer wrapper classes by flattening the
286+
state dicts from each object into one state dict before saving/loading.
287+
We rely on the individual state_dicts to not collide, which is gauranteed for
288+
the model by correct pipeline splitting and for the optimizer by the flattening
289+
support described in (1).
290+
291+
3. LR schedulers also index model states like optimizers and would need to be
292+
flattened properly to support resharding. Unfortunately, the implementations of
293+
different lr_schedulers do not follow a clear pattern like optimizers do, so it's
294+
hard to write a generic 'flattener' utility.
295+
296+
TODO: This is currently unsolved and needs a fix.
297+
"""
298+
self.states = states
299+
self.states.update(
300+
{
301+
"model": ModelWrapper(model_parts),
302+
"optimizer": optimizers,
303+
"dataloader": dataloader,
304+
"lr_scheduler": lr_schedulers,
305+
}
306+
)
307+
245308
def _create_checkpoint_id(self, step: int) -> str:
246309
return os.path.join(self.folder, f"step-{step}")
247310

@@ -324,31 +387,8 @@ def _async_wait(self) -> None:
324387
self.async_future.result()
325388

326389
def _async_with_pinned_memory(self, checkpoint_id: str) -> None:
327-
try:
328-
from torch.distributed._state_dict_utils import (
329-
_copy_state_dict,
330-
_create_cpu_state_dict,
331-
)
332-
except ImportError as e:
333-
raise ImportError(
334-
"Please install the latest PyTorch nightly to use async checkpointing with pinned memory."
335-
) from e
336-
state_dict = dcp.state_dict_saver._stateful_to_state_dict(self.states)
337-
if self.cpu_offload_state_dict is None:
338-
logger.debug(f"Preparing the CPU memory, {time.monotonic()=}.:.2f")
339-
self.cpu_offload_state_dict = _create_cpu_state_dict(
340-
state_dict, pin_memory=True, share_memory=True
341-
)
342-
343-
logger.debug(f"Staging the state_dict, {time.monotonic()=}.:.2f")
344-
with torch.cuda.stream(self.staging_stream):
345-
self.cpu_offload_state_dict = _copy_state_dict(
346-
state_dict,
347-
self.cpu_offload_state_dict,
348-
non_blocking=True,
349-
)
350-
self.staging = True
351-
self.staging_id = checkpoint_id
390+
self.cpu_staging(checkpoint_id)
391+
self.sending_to_checkpoint_mp = True
352392

353393
def save(self, curr_step: int, force: bool = False) -> None:
354394
"""
@@ -358,6 +398,8 @@ def save(self, curr_step: int, force: bool = False) -> None:
358398
for initial seed checkpoint.
359399
"""
360400
if not self._should_save(curr_step, force):
401+
if self.ft_manager:
402+
self.cpu_staging(None)
361403
return
362404

363405
begin = time.monotonic()
@@ -381,26 +423,51 @@ def save(self, curr_step: int, force: bool = False) -> None:
381423
f"in {time.monotonic() - begin:.2f} seconds."
382424
)
383425

426+
def cpu_staging(self, checkpoint_id: Optional[str]) -> None:
427+
"""Offload state_dict to CPU memory"""
428+
state_dict = dcp.state_dict_saver._stateful_to_state_dict(self.states)
429+
if self.cpu_offload_state_dict is None:
430+
logger.debug(f"Preparing the CPU memory, {time.monotonic()=}.:.2f")
431+
self.cpu_offload_state_dict = _create_cpu_state_dict(
432+
state_dict, pin_memory=True, share_memory=True
433+
)
434+
435+
logger.debug(f"Staging the state_dict, {time.monotonic()=}.:.2f")
436+
with torch.cuda.stream(self.staging_stream):
437+
self.cpu_offload_state_dict = _copy_state_dict(
438+
state_dict,
439+
self.cpu_offload_state_dict,
440+
non_blocking=True,
441+
)
442+
self.staging = True
443+
self.staging_id = checkpoint_id
444+
445+
def wait_for_staging(self) -> None:
446+
if not self.staging_stream.query():
447+
self.staging_stream.synchronize()
448+
self.staging = False
449+
450+
def staging_results(self) -> Dict[str, Any]:
451+
self.maybe_wait_for_staging()
452+
return self.cpu_offload_state_dict
453+
384454
def maybe_wait_for_staging(self) -> None:
385-
if (
386-
self.enable_checkpoint
387-
and self.async_mode == AsyncMode.ASYNC_WITH_PINNED_MEM
388-
and self.staging
389-
):
390-
if not self.staging_stream.query():
391-
self.staging_stream.synchronize()
392-
393-
def sync_func():
394-
self.mp_queue_send.put_nowait(
395-
(self.cpu_offload_state_dict, self.staging_id)
396-
)
397-
398-
# This may be a faster way to do zero-overhead checkpointing staging
399-
# checkpointing but we need more thorough investigation before
400-
# swithing to this method.
401-
# self.my_thread = threading.Thread(target=func).start()
402-
sync_func()
403-
self.staging = False
455+
if self.enable_staging and self.staging:
456+
self.wait_for_staging()
457+
458+
if self.sending_to_checkpoint_mp:
459+
# Copy the sync staging result to another process.
460+
def sync_func():
461+
self.mp_queue_send.put_nowait(
462+
(self.cpu_offload_state_dict, self.staging_id)
463+
)
464+
465+
# This may be a faster way to do zero-overhead checkpointing staging
466+
# checkpointing but we need more thorough investigation before
467+
# swithing to this method.
468+
# self.my_thread = threading.Thread(target=func).start()
469+
sync_func()
470+
self.sending_to_checkpoint_mp = False
404471

405472
def load(self, step: int = -1) -> bool:
406473
if not self.enable_checkpoint:

‎torchtitan/config_manager.py

+13
Original file line numberDiff line numberDiff line change
@@ -631,6 +631,19 @@ def __init__(self):
631631
action="store_true",
632632
)
633633

634+
self.parser.add_argument(
635+
"--experimental.enable_torchft",
636+
action="store_true",
637+
help="Enable TorchFT integration.",
638+
)
639+
640+
self.parser.add_argument(
641+
"--experimental.ft_replica_group_id",
642+
type=int,
643+
default=-1,
644+
help="The FT replicate group of this run.",
645+
)
646+
634647
def to_dict(self):
635648
return self.args_dict
636649

‎torchtitan/ft.py

+58
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,58 @@
1+
import importlib
2+
from typing import Any, Callable, Optional
3+
4+
from torch.distributed._state_dict_utils import _copy_state_dict, _create_cpu_state_dict
5+
6+
from torchtitan.config_manager import JobConfig
7+
8+
if importlib.util.find_spec("torchft") is not None:
9+
import torchft as ft
10+
11+
has_torchft = True
12+
else:
13+
has_torchft = False
14+
15+
16+
def init_ft_manager(job: JobConfig) -> Optional["ft.Manager"]:
17+
"""
18+
Initialize the FT manager for the given job.
19+
"""
20+
if not job.experimental.enable_torchft:
21+
return None
22+
23+
if not has_torchft:
24+
raise ImportError("torchft is not installed. Please install it.")
25+
26+
pg = ft.ProcessGroupBabyNCCL()
27+
manager = ft.Manager(
28+
pg=pg,
29+
min_replica_size=1,
30+
load_state_dict=None,
31+
state_dict=None,
32+
use_async_quorum=True,
33+
replica_id=f"torchtitan_ft_{job.experimental.ft_replica_group_id}",
34+
)
35+
36+
return manager
37+
38+
39+
def set_ft_state_dict_fns(manager: Optional["ft.Manager"], ckpt_manager) -> None:
40+
"""
41+
Set the state dict for the given manager.
42+
"""
43+
if manager is None:
44+
return
45+
46+
def state_dict():
47+
ret = {}
48+
for k, v in ckpt_manager.staging_results().items():
49+
if k in {"model", "optimizer", "lr_schedulers"}:
50+
ret[k] = v
51+
return ret
52+
53+
def load_state_dict(state_dict):
54+
assert state_dict is not None
55+
for k, v in state_dict.items():
56+
ckpt_manager.states[k].load_state_dict(v)
57+
58+
manager.set_state_dict_fns(load_state_dict, state_dict)

‎torchtitan/optimizer.py

+33-6
Original file line numberDiff line numberDiff line change
@@ -177,8 +177,32 @@ def zero_grad(self) -> None:
177177
pass
178178

179179

180+
class FTOptimizersContainer(Optimizer):
181+
def __init__(
182+
self,
183+
model_parts: List[nn.Module],
184+
optimizer_kwargs: Dict[str, Any],
185+
name: str,
186+
ft_manager: Any,
187+
) -> None:
188+
import torchft as ft
189+
190+
super().__init__()
191+
192+
# Force to initialize the optimizer state so that `optim.step()`
193+
# won't be called by state_dict() and load_state_dict().
194+
_ = {
195+
k: v
196+
for sd in map(get_optimizer_state_dict, model_parts, self.optimizers)
197+
for k, v in sd.items()
198+
}
199+
self.optimizers = [ft.Optimizer(ft_manager, optim) for optim in self.optimizers]
200+
201+
180202
def build_optimizers(
181-
model_parts: List[nn.Module], job_config: JobConfig
203+
model_parts: List[nn.Module],
204+
job_config: JobConfig,
205+
ft_manager: Optional[Any] = None,
182206
) -> OptimizersContainer:
183207
"""Create a OptimizersContainer for the given model parts and job config.
184208
@@ -213,11 +237,14 @@ def build_optimizers(
213237
"foreach": not fused,
214238
}
215239

216-
return (
217-
OptimizersContainer(model_parts, optimizer_kwargs, name)
218-
if not optim_in_bwd
219-
else OptimizersInBackwardContainer(model_parts, optimizer_kwargs, name)
220-
)
240+
if optim_in_bwd and ft_manager:
241+
raise ValueError("TorchFT is not supported with optimizers in backward.")
242+
elif optim_in_bwd:
243+
return OptimizersInBackwardContainer(model_parts, optimizer_kwargs, name)
244+
elif ft_manager:
245+
return FTOptimizersContainer(model_parts, optimizer_kwargs, name, ft_manager)
246+
else:
247+
return OptimizersContainer(model_parts, optimizer_kwargs, name)
221248

222249

223250
class LRSchedulersContainer(Stateful):

0 commit comments

Comments
 (0)
Failed to load comments.