communication_op.py 7.16 KB
Newer Older
1
2
3
from collections import namedtuple
from typing import Any, Dict, List, Optional, Union

4
5
from torch.distributed import ProcessGroup

6
7
8
import torch

from vllm.model_executor.parallel_utils.parallel_state import (
9
    get_tensor_model_parallel_rank,
10
11
12
13
14
    get_tensor_model_parallel_world_size,
    get_tensor_model_parallel_group,
)


15
def tensor_model_parallel_all_reduce(input_: torch.Tensor) -> torch.Tensor:
16
17
    """All-reduce the input tensor across model parallel group.

18
    NOTE: This operation is applied in-place on the input tensor.
19
20
21
22
23
    """
    # Bypass the function if we are using only 1 GPU.
    if get_tensor_model_parallel_world_size() == 1:
        return input_
    # All-reduce.
24
25
    torch.distributed.all_reduce(input_,
                                 group=get_tensor_model_parallel_group())
26
27
28
    return input_


29
30
def tensor_model_parallel_all_gather(input_: torch.Tensor,
                                     dim: int = -1) -> torch.Tensor:
31
32
33
34
35
36
37
38
39
40
41
42
43
44
45
46
47
48
49
50
51
52
53
54
    """All-gather the input tensor across model parallel group."""
    world_size = get_tensor_model_parallel_world_size()
    # Bypass the function if we are using only 1 GPU.
    if world_size == 1:
        return input_
    assert -input_.dim() <= dim < input_.dim(), (
        f"Invalid dim ({dim}) for input tensor with shape {input_.size()}")
    if dim < 0:
        # Convert negative dim to positive.
        dim += input_.dim()
    input_size = input_.size()
    # Allocate output tensor.
    output_tensor = torch.empty((world_size, ) + input_size,
                                dtype=input_.dtype,
                                device=input_.device)
    # All-gather.
    torch.distributed.all_gather_into_tensor(
        output_tensor, input_, group=get_tensor_model_parallel_group())
    # Reshape
    output_tensor = output_tensor.movedim(0, dim)
    output_tensor = output_tensor.reshape(input_size[:dim] +
                                          (world_size * input_size[dim], ) +
                                          input_size[dim + 1:])
    return output_tensor
55
56


57
58
59
def tensor_model_parallel_gather(input_: torch.Tensor,
                                 dst: int = 0,
                                 dim: int = -1) -> torch.Tensor:
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
    """Gather the input tensor across model parallel group.

    NOTE: We assume that the input tensor is on the same device across
    all the ranks.
    """
    world_size = get_tensor_model_parallel_world_size()
    # Bypass the function if we are using only 1 GPU.
    if world_size == 1:
        return input_
    assert -input_.dim() <= dim < input_.dim(), (
        f"Invalid dim ({dim}) for input tensor with shape {input_.size()}")
    if dim < 0:
        # Convert negative dim to positive.
        dim += input_.dim()
    # Allocate output tensor.
    if get_tensor_model_parallel_rank() == dst:
        gather_list = [torch.empty_like(input_) for _ in range(world_size)]
    else:
        gather_list = None
    # Gather.
    torch.distributed.gather(input_,
                             gather_list,
                             dst=dst,
                             group=get_tensor_model_parallel_group())
    if get_tensor_model_parallel_rank() == dst:
        output_tensor = torch.cat(gather_list, dim=dim)
    else:
        output_tensor = None
    return output_tensor


91
92
93
def broadcast(input_: torch.Tensor,
              src: int = 0,
              group: Optional[ProcessGroup] = None):
94
    """Broadcast the input tensor."""
95
96
97
    group = group or torch.distributed.group.WORLD
    ranks = torch.distributed.get_process_group_ranks(group)
    assert src in ranks, f"Invalid src rank ({src})"
98
99

    # Bypass the function if we are using only 1 GPU.
100
    world_size = torch.distributed.get_world_size(group=group)
101
102
103
    if world_size == 1:
        return input_
    # Broadcast.
104
    torch.distributed.broadcast(input_, src=src, group=group)
105
106
107
    return input_


108
109
110
def broadcast_object_list(obj_list: List[Any],
                          src: int = 0,
                          group: Optional[ProcessGroup] = None):
111
    """Broadcast the input object list."""
112
113
114
    group = group or torch.distributed.group.WORLD
    ranks = torch.distributed.get_process_group_ranks(group)
    assert src in ranks, f"Invalid src rank ({src})"
115
116

    # Bypass the function if we are using only 1 GPU.
117
    world_size = torch.distributed.get_world_size(group=group)
118
119
120
    if world_size == 1:
        return obj_list
    # Broadcast.
121
    torch.distributed.broadcast_object_list(obj_list, src=src, group=group)
122
    return obj_list
123
124
125
126
127


TensorMetadata = namedtuple("TensorMetadata", ["dtype", "size"])


128
129
130
131
132
def broadcast_tensor_dict(
    tensor_dict: Optional[Dict[Any, Union[torch.Tensor, Any]]] = None,
    src: int = 0,
    group: Optional[ProcessGroup] = None,
) -> Dict[Any, Union[torch.Tensor, Any]]:
133
    """Broadcast the input tensor dictionary."""
134
135
136
    group = group or torch.distributed.group.WORLD
    ranks = torch.distributed.get_process_group_ranks(group)
    assert src in ranks, f"Invalid src rank ({src})"
137
138

    # Bypass the function if we are using only 1 GPU.
139
    world_size = torch.distributed.get_world_size(group=group)
140
141
142
    if world_size == 1:
        return tensor_dict

143
    rank = torch.distributed.get_rank()
144
145
146
147
148
149
150
151
152
153
154
155
156
157
    if rank == src:
        assert isinstance(
            tensor_dict,
            dict), (f"Expecting a dictionary, got {type(tensor_dict)}")
        metadata_list = []
        for key, value in tensor_dict.items():
            if isinstance(value, torch.Tensor):
                assert value.is_cuda, (
                    f"Tensor {key}: {value} is not on cuda. Currently we only "
                    f"support broadcasting tensors on cuda.")
                metadata_list.append(
                    (key, TensorMetadata(value.dtype, value.size())))
            else:
                metadata_list.append((key, value))
158
159
160
        torch.distributed.broadcast_object_list([metadata_list],
                                                src=src,
                                                group=group)
161
162
163
164
165
166
        for key, value in metadata_list:
            if isinstance(value, TensorMetadata):
                tensor = tensor_dict[key]
                torch.distributed.broadcast(tensor, src=src)
    else:
        recv_metadata_list = [None]
167
168
169
        torch.distributed.broadcast_object_list(recv_metadata_list,
                                                src=src,
                                                group=group)
170
171
172
173
174
175
176
177
178
179
        metadata_list = recv_metadata_list[0]
        tensor_dict = {}
        async_handles = []
        for key, value in metadata_list:
            if isinstance(value, TensorMetadata):
                tensor = torch.empty(value.size,
                                     dtype=value.dtype,
                                     device="cuda")
                async_handle = torch.distributed.broadcast(tensor,
                                                           src=src,
180
181
                                                           async_op=True,
                                                           group=group)
182
183
184
185
186
187
188
                async_handles.append(async_handle)
                tensor_dict[key] = tensor
            else:
                tensor_dict[key] = value
        for async_handle in async_handles:
            async_handle.wait()
    return tensor_dict