mirror of
https://github.com/zebrajr/pytorch.git
synced 2025-12-07 12:21:27 +01:00
Summary: Pull Request resolved: https://github.com/pytorch/pytorch/pull/40222 Mention the TensorPipe agent in the RPC docs and give users the information they need to choose which agent to use. ghstack-source-id: 106225711 Test Plan: Export to GitHub, build locally and try out the docs. Differential Revision: D22116494 fbshipit-source-id: 30703ba8410c40f64e785f60d71dfd9faa8de4a1
207 lines
6.3 KiB
Python
207 lines
6.3 KiB
Python
from __future__ import absolute_import, division, print_function, unicode_literals
|
|
|
|
import collections
|
|
from datetime import timedelta
|
|
import enum
|
|
|
|
import torch.distributed as dist
|
|
|
|
from . import constants as rpc_constants
|
|
|
|
|
|
BackendValue = collections.namedtuple(
|
|
"BackendValue", ["construct_rpc_backend_options_handler", "init_backend_handler"]
|
|
)
|
|
|
|
|
|
def _backend_type_repr(self):
|
|
return "BackendType." + self.name
|
|
|
|
|
|
_backend_type_doc = """
|
|
An enum class of available backends.
|
|
|
|
PyTorch ships with two builtin backends: ``BackendType.PROCESS_GROUP`` and
|
|
``BackendType.TENSORPIPE``. Additional ones can be registered using the
|
|
:func:`~torch.distributed.rpc.backend_registry.register_backend` function.
|
|
"""
|
|
|
|
# Create an enum type, `BackendType`, with empty members.
|
|
BackendType = enum.Enum(value="BackendType", names={})
|
|
BackendType.__repr__ = _backend_type_repr
|
|
BackendType.__doc__ = _backend_type_doc
|
|
|
|
def backend_registered(backend_name):
|
|
"""
|
|
Checks if backend_name is registered as an RPC backend.
|
|
|
|
Arguments:
|
|
backend_name (str): string to identify the RPC backend.
|
|
Returns:
|
|
True if the backend has been registered with ``register_backend``, else
|
|
False.
|
|
"""
|
|
return backend_name in BackendType.__members__.keys()
|
|
|
|
|
|
def register_backend(
|
|
backend_name, construct_rpc_backend_options_handler, init_backend_handler
|
|
):
|
|
"""Registers a new RPC backend.
|
|
|
|
Arguments:
|
|
backend_name (str): backend string to identify the handler.
|
|
construct_rpc_backend_options_handler (function):
|
|
Handler that is invoked when
|
|
rpc_backend.construct_rpc_backend_options(**dict) is called.
|
|
init_backend_handler (function): Handler that is invoked when the
|
|
`_init_rpc_backend()` function is called with a backend.
|
|
This returns the agent.
|
|
"""
|
|
global BackendType
|
|
if backend_registered(backend_name):
|
|
raise RuntimeError("RPC backend {}: already registered".format(backend_name))
|
|
# Create a new enum type, `BackendType`, with extended members.
|
|
existing_enum_dict = {member.name: member.value for member in BackendType}
|
|
extended_enum_dict = dict(
|
|
{
|
|
backend_name: BackendValue(
|
|
construct_rpc_backend_options_handler=construct_rpc_backend_options_handler,
|
|
init_backend_handler=init_backend_handler,
|
|
)
|
|
},
|
|
**existing_enum_dict
|
|
)
|
|
BackendType = enum.Enum(value="BackendType", names=extended_enum_dict)
|
|
BackendType.__repr__ = _backend_type_repr
|
|
BackendType.__doc__ = _backend_type_doc
|
|
return BackendType[backend_name]
|
|
|
|
|
|
def construct_rpc_backend_options(
|
|
backend,
|
|
rpc_timeout=rpc_constants.DEFAULT_RPC_TIMEOUT_SEC,
|
|
init_method=rpc_constants.DEFAULT_INIT_METHOD,
|
|
**kwargs
|
|
):
|
|
|
|
return backend.value.construct_rpc_backend_options_handler(
|
|
rpc_timeout, init_method, **kwargs
|
|
)
|
|
|
|
|
|
def init_backend(backend, *args, **kwargs):
|
|
return backend.value.init_backend_handler(*args, **kwargs)
|
|
|
|
|
|
def _process_group_construct_rpc_backend_options_handler(
|
|
rpc_timeout,
|
|
init_method,
|
|
num_send_recv_threads=rpc_constants.DEFAULT_NUM_SEND_RECV_THREADS,
|
|
**kwargs
|
|
):
|
|
from . import ProcessGroupRpcBackendOptions
|
|
|
|
return ProcessGroupRpcBackendOptions(
|
|
rpc_timeout=rpc_timeout,
|
|
init_method=init_method,
|
|
num_send_recv_threads=num_send_recv_threads
|
|
)
|
|
|
|
def _init_process_group(store, rank, world_size):
|
|
# Initialize ProcessGroup.
|
|
process_group_timeout = rpc_constants.DEFAULT_PROCESS_GROUP_TIMEOUT
|
|
|
|
# We're using a bunch of private APIs here since `new_group` requires the
|
|
# default group to be initialized.
|
|
group = dist.ProcessGroupGloo(store, rank, world_size, process_group_timeout)
|
|
|
|
assert group is not None, "Failed to initialize default ProcessGroup."
|
|
|
|
if (rank != -1) and (rank != group.rank()):
|
|
raise RuntimeError(
|
|
"rank argument {} doesn't match pg rank {}".format(rank, group.rank())
|
|
)
|
|
if (world_size != -1) and (world_size != group.size()):
|
|
raise RuntimeError(
|
|
"world_size argument {} doesn't match pg size {}".format(
|
|
world_size, group.size()
|
|
)
|
|
)
|
|
return group
|
|
|
|
def _process_group_init_backend_handler(
|
|
store, name, rank, world_size, rpc_backend_options
|
|
):
|
|
from . import ProcessGroupAgent
|
|
|
|
group = _init_process_group(store, rank, world_size)
|
|
|
|
# TODO: add try-except and destroy _agent in all processes if any fails.
|
|
return ProcessGroupAgent(
|
|
name,
|
|
group,
|
|
rpc_backend_options.num_send_recv_threads,
|
|
timedelta(seconds=rpc_backend_options.rpc_timeout),
|
|
)
|
|
|
|
|
|
register_backend(
|
|
"PROCESS_GROUP",
|
|
_process_group_construct_rpc_backend_options_handler,
|
|
_process_group_init_backend_handler,
|
|
)
|
|
|
|
def _tensorpipe_construct_rpc_backend_options_handler(
|
|
rpc_timeout,
|
|
init_method,
|
|
num_worker_threads=rpc_constants.DEFAULT_NUM_WORKER_THREADS,
|
|
_transports=None,
|
|
_channels=None,
|
|
**kwargs
|
|
):
|
|
from . import TensorPipeRpcBackendOptions
|
|
|
|
return TensorPipeRpcBackendOptions(
|
|
rpc_timeout=rpc_timeout,
|
|
init_method=init_method,
|
|
num_worker_threads=num_worker_threads,
|
|
_transports=_transports,
|
|
_channels=_channels,
|
|
)
|
|
|
|
|
|
def _tensorpipe_init_backend_handler(store, name, rank, world_size, rpc_backend_options):
|
|
from . import TensorPipeRpcBackendOptions
|
|
from . import TensorPipeAgent
|
|
|
|
if not isinstance(store, dist.Store):
|
|
raise TypeError("`store` must be a c10d::Store. {}".format(store))
|
|
|
|
if not isinstance(
|
|
rpc_backend_options, TensorPipeRpcBackendOptions
|
|
):
|
|
raise TypeError(
|
|
"`rpc_backend_options` must be a `TensorPipeRpcBackendOptions`. {}".format(
|
|
rpc_backend_options
|
|
)
|
|
)
|
|
|
|
# The agent's join method is required to behave like a barrier and perform
|
|
# collective operations, for which it relies on a process group, instead of
|
|
# re-implementing this on top of RPCs.
|
|
|
|
group = _init_process_group(store, rank, world_size)
|
|
|
|
# TODO: add try-except and destroy _agent in all processes if any fails.
|
|
return TensorPipeAgent(
|
|
store, name, rank, world_size, group, rpc_backend_options
|
|
)
|
|
|
|
|
|
register_backend(
|
|
"TENSORPIPE",
|
|
_tensorpipe_construct_rpc_backend_options_handler,
|
|
_tensorpipe_init_backend_handler,
|
|
)
|