• 文档 >
  • 循环DQN:训练循环策略
快捷方式

循环DQN:训练循环策略

**作者**:Vincent Moens

您将学到什么
  • 如何在 TorchRL 中将 RNN 整合到 Actor 中

  • 如何将基于内存的策略与回放缓冲区和损失模块一起使用

先决条件
  • PyTorch v2.0.0

  • gym[mujoco]

  • tqdm

概述

基于内存的策略不仅在观测值部分可观测时至关重要,而且在必须考虑时间维度才能做出明智决策时也至关重要。

循环神经网络长期以来一直是基于内存的策略的常用工具。其思想是在两个连续步骤之间将循环状态保存在内存中,并将其与当前观测值一起用作策略的输入。

本教程展示了如何在使用 TorchRL 的策略中整合 RNN。

主要学习内容

  • 在 TorchRL 中将 RNN 整合到 Actor 中;

  • 将基于内存的策略与回放缓冲区和损失模块一起使用。

在 TorchRL 中使用 RNN 的核心思想是使用 TensorDict 作为数据载体,将隐藏状态从一步传递到下一步。我们将构建一个策略,该策略从当前 TensorDict 中读取先前的循环状态,并在下一个状态的 TensorDict 中写入当前循环状态

Data collection with a recurrent policy

如图所示,我们的环境使用归零的循环状态填充 TensorDict,这些状态与观测值一起被策略读取以生成动作,以及将在下一步中使用的循环状态。当调用step_mdp() 函数时,下一个状态的循环状态将被带到当前 TensorDict。让我们看看在实践中是如何实现的。

如果您在 Google Colab 中运行此代码,请确保安装以下依赖项

!pip3 install torchrl
!pip3 install gym[mujoco]
!pip3 install tqdm

设置

import torch
import tqdm
from tensordict.nn import (
    TensorDictModule as Mod,
    TensorDictSequential,
    TensorDictSequential as Seq,
)
from torch import nn
from torchrl.collectors import SyncDataCollector
from torchrl.data import LazyMemmapStorage, TensorDictReplayBuffer
from torchrl.envs import (
    Compose,
    ExplorationType,
    GrayScale,
    InitTracker,
    ObservationNorm,
    Resize,
    RewardScaling,
    set_exploration_type,
    StepCounter,
    ToTensorImage,
    TransformedEnv,
)
from torchrl.envs.libs.gym import GymEnv
from torchrl.modules import ConvNet, EGreedyModule, LSTMModule, MLP, QValueModule
from torchrl.objectives import DQNLoss, SoftUpdate

is_fork = multiprocessing.get_start_method() == "fork"
device = (
    torch.device(0)
    if torch.cuda.is_available() and not is_fork
    else torch.device("cpu")
)

环境

像往常一样,第一步是构建我们的环境:它帮助我们定义问题并相应地构建策略网络。在本教程中,我们将运行 CartPole gym 环境的单个基于像素的实例,并使用一些自定义转换:转换为灰度、调整大小到 84x84、缩小奖励并标准化观测值。

注意

StepCounter 转换是辅助性的。由于 CartPole 任务的目标是使轨迹尽可能长,因此计算步数可以帮助我们跟踪策略的性能。

两个转换对于本教程的目的很重要

  • InitTracker 将通过在 TensorDict 中添加一个"is_init" 布尔掩码来标记对reset() 的调用,这将跟踪哪些步骤需要重置 RNN 隐藏状态。

  • TensorDictPrimer 转换稍微复杂一些。它不是使用 RNN 策略所必需的。但是,它指示环境(以及随后的收集器)应该预期一些额外的键。添加后,对env.reset() 的调用将使用归零的张量填充 primer 中指示的条目。知道策略期望这些张量,收集器将在收集期间将其传递。最终,我们将循环状态存储在回放缓冲区中,这将帮助我们引导损失模块中 RNN 操作的计算(否则将以 0s 初始化)。总之:不包含此转换不会对策略的训练产生重大影响,但它将导致收集的数据和回放缓冲区中的循环键消失,这反过来又会导致训练效果略微不佳。幸运的是,我们提出的LSTMModule 配备了一个辅助方法来为我们构建转换,因此我们可以等到构建它!

env = TransformedEnv(
    GymEnv("CartPole-v1", from_pixels=True, device=device),
    Compose(
        ToTensorImage(),
        GrayScale(),
        Resize(84, 84),
        StepCounter(),
        InitTracker(),
        RewardScaling(loc=0.0, scale=0.1),
        ObservationNorm(standard_normal=True, in_keys=["pixels"]),
    ),
)

像往常一样,我们需要手动初始化我们的标准化常数

env.transform[-1].init_stats(1000, reduce_dim=[0, 1, 2], cat_dim=0, keep_dims=[0])
td = env.reset()

策略

我们的策略将包含三个组件:ConvNet 主干、LSTMModule 内存层和一个浅层MLP 块,它将 LSTM 输出映射到动作值。

卷积神经网络

我们构建了一个卷积神经网络,两侧带有torch.nn.AdaptiveAvgPool2d,它将输出压缩到大小为 64 的向量中。ConvNet 可以帮助我们做到这一点

feature = Mod(
    ConvNet(
        num_cells=[32, 32, 64],
        squeeze_output=True,
        aggregator_class=nn.AdaptiveAvgPool2d,
        aggregator_kwargs={"output_size": (1, 1)},
        device=device,
    ),
    in_keys=["pixels"],
    out_keys=["embed"],
)

我们在数据批次上执行第一个模块以收集输出向量的尺寸

n_cells = feature(env.reset())["embed"].shape[-1]

LSTM 模块

TorchRL 提供了一个专门的LSTMModule 类来将 LSTM 整合到您的代码库中。它是一个TensorDictModuleBase 子类:因此,它有一组in_keysout_keys,指示在模块执行期间应读取和写入/更新哪些值。该类附带可自定义的预定义值,以方便其构建。

注意

使用限制:该类支持几乎所有 LSTM 功能,例如 dropout 或多层 LSTM。但是,为了遵守 TorchRL 的约定,此 LSTM 必须将batch_first 属性设置为True,这**不是** PyTorch 中的默认值。但是,我们的LSTMModule 更改了此默认行为,因此我们可以使用原生调用。

此外,LSTM 不能将bidirectional 属性设置为True,因为这在联机设置中不可用。在这种情况下,默认值是正确的。

lstm = LSTMModule(
    input_size=n_cells,
    hidden_size=128,
    device=device,
    in_key="embed",
    out_key="embed",
)

让我们看一下 LSTM 模块类,特别是它的 in 和 out_keys

print("in_keys", lstm.in_keys)
print("out_keys", lstm.out_keys)
in_keys ['embed', 'recurrent_state_h', 'recurrent_state_c', 'is_init']
out_keys ['embed', ('next', 'recurrent_state_h'), ('next', 'recurrent_state_c')]

我们可以看到这些值包含我们指示为 in_key(和 out_key)的键以及循环键名称。out_keys 前面带有“next”前缀,表示它们需要写入“next”TensorDict 中。我们使用此约定(可以通过传递 in_keys/out_keys 参数来覆盖),以确保对 step_mdp() 的调用会将循环状态移动到根 TensorDict,使其在后续调用期间可供 RNN 使用(请参阅引言中的图)。

如前所述,我们还有一个可选的转换要添加到我们的环境中,以确保将循环状态传递到缓冲区。 make_tensordict_primer() 方法正是这样做的

env.append_transform(lstm.make_tensordict_primer())
TransformedEnv(
    env=GymEnv(env=CartPole-v1, batch_size=torch.Size([]), device=cpu),
    transform=Compose(
            ToTensorImage(keys=['pixels']),
            GrayScale(keys=['pixels']),
            Resize(w=84, h=84, interpolation=InterpolationMode.BILINEAR, keys=['pixels']),
            StepCounter(keys=[]),
            InitTracker(keys=[]),
            RewardScaling(loc=0.0000, scale=0.1000, keys=['reward']),
            ObservationNorm(keys=['pixels']),
            TensorDictPrimer(primers=CompositeSpec(
                recurrent_state_h: UnboundedContinuousTensorSpec(
                    shape=torch.Size([1, 128]),
                    space=None,
                    device=cpu,
                    dtype=torch.float32,
                    domain=continuous),
                recurrent_state_c: UnboundedContinuousTensorSpec(
                    shape=torch.Size([1, 128]),
                    space=None,
                    device=cpu,
                    dtype=torch.float32,
                    domain=continuous),
                device=cpu,
                shape=torch.Size([])), default_value={'recurrent_state_h': 0.0, 'recurrent_state_c': 0.0}, random=None)))

就是这样!我们可以在添加了 primer 后打印环境以检查一切是否正常。

print(env)
TransformedEnv(
    env=GymEnv(env=CartPole-v1, batch_size=torch.Size([]), device=cpu),
    transform=Compose(
            ToTensorImage(keys=['pixels']),
            GrayScale(keys=['pixels']),
            Resize(w=84, h=84, interpolation=InterpolationMode.BILINEAR, keys=['pixels']),
            StepCounter(keys=[]),
            InitTracker(keys=[]),
            RewardScaling(loc=0.0000, scale=0.1000, keys=['reward']),
            ObservationNorm(keys=['pixels']),
            TensorDictPrimer(primers=CompositeSpec(
                recurrent_state_h: UnboundedContinuousTensorSpec(
                    shape=torch.Size([1, 128]),
                    space=None,
                    device=cpu,
                    dtype=torch.float32,
                    domain=continuous),
                recurrent_state_c: UnboundedContinuousTensorSpec(
                    shape=torch.Size([1, 128]),
                    space=None,
                    device=cpu,
                    dtype=torch.float32,
                    domain=continuous),
                device=cpu,
                shape=torch.Size([])), default_value={'recurrent_state_h': 0.0, 'recurrent_state_c': 0.0}, random=None)))

MLP

我们使用单层 MLP 来表示我们将用于策略的动作值。

mlp = MLP(
    out_features=2,
    num_cells=[
        64,
    ],
    device=device,
)

并用零填充偏差。

mlp[-1].bias.data.fill_(0.0)
mlp = Mod(mlp, in_keys=["embed"], out_keys=["action_value"])

使用 Q 值选择动作

我们策略的最后一部分是 Q 值模块。Q 值模块 QValueModule 将读取我们的 MLP 生成的 "action_values" 键,并从中收集具有最大值的动作。我们唯一需要做的是指定动作空间,这可以通过传递字符串或动作规范来完成。这使我们能够使用类别(有时称为“稀疏”)编码或其单热版本。

qval = QValueModule(action_space=None, spec=env.action_spec)

注意

TorchRL 还提供了一个包装类 torchrl.modules.QValueActor,它将一个模块与 QValueModule 一起包装在一个 Sequential 中,就像我们在这里显式地做的那样。这样做几乎没有优势,而且过程也不够透明,但最终结果将与我们在这里所做的类似。

我们现在可以将它们放在 TensorDictSequential 中。

stoch_policy = Seq(feature, lstm, mlp, qval)

DQN 是一种确定性算法,探索是其至关重要的部分。我们将使用 \(\epsilon\)-贪婪策略,其中 epsilon 为 0.2,逐渐衰减到 0。这种衰减是通过调用 step() 实现的(请参阅下面的训练循环)。

exploration_module = EGreedyModule(
    annealing_num_steps=1_000_000, spec=env.action_spec, eps_init=0.2
)
stoch_policy = TensorDictSequential(
    stoch_policy,
    exploration_module,
)

使用模型进行损失

我们构建的模型非常适合在顺序设置中使用。但是,类 torch.nn.LSTM 可以使用 cuDNN 优化的后端在 GPU 设备上更快地运行 RNN 序列。我们不希望错过这样的加速训练循环的机会!要使用它,我们只需要告诉 LSTM 模块在损失使用时以“循环模式”运行。由于我们通常希望拥有 LSTM 模块的两个副本,因此我们通过调用 set_recurrent_mode() 方法来实现,该方法将返回 LSTM 的一个新实例(具有共享权重),该实例将假设输入数据本质上是顺序的。

policy = Seq(feature, lstm.set_recurrent_mode(True), mlp, qval)

因为我们仍然有一些未初始化的参数,所以我们应该在创建优化器等之前初始化它们。

policy(env.reset())
TensorDict(
    fields={
        action: Tensor(shape=torch.Size([2]), device=cpu, dtype=torch.int64, is_shared=False),
        action_value: Tensor(shape=torch.Size([2]), device=cpu, dtype=torch.float32, is_shared=False),
        chosen_action_value: Tensor(shape=torch.Size([1]), device=cpu, dtype=torch.float32, is_shared=False),
        done: Tensor(shape=torch.Size([1]), device=cpu, dtype=torch.bool, is_shared=False),
        embed: Tensor(shape=torch.Size([128]), device=cpu, dtype=torch.float32, is_shared=False),
        is_init: Tensor(shape=torch.Size([1]), device=cpu, dtype=torch.bool, is_shared=False),
        next: TensorDict(
            fields={
                recurrent_state_c: Tensor(shape=torch.Size([1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
                recurrent_state_h: Tensor(shape=torch.Size([1, 128]), device=cpu, dtype=torch.float32, is_shared=False)},
            batch_size=torch.Size([]),
            device=cpu,
            is_shared=False),
        pixels: Tensor(shape=torch.Size([1, 84, 84]), device=cpu, dtype=torch.float32, is_shared=False),
        recurrent_state_c: Tensor(shape=torch.Size([1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
        recurrent_state_h: Tensor(shape=torch.Size([1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
        step_count: Tensor(shape=torch.Size([1]), device=cpu, dtype=torch.int64, is_shared=False),
        terminated: Tensor(shape=torch.Size([1]), device=cpu, dtype=torch.bool, is_shared=False),
        truncated: Tensor(shape=torch.Size([1]), device=cpu, dtype=torch.bool, is_shared=False)},
    batch_size=torch.Size([]),
    device=cpu,
    is_shared=False)

DQN 损失

我们的 DQN 损失要求我们传递策略,以及动作空间。虽然这看起来可能是多余的,但它很重要,因为我们希望确保 DQNLossQValueModule 类兼容,但彼此之间没有强依赖关系。

要使用 Double-DQN,我们要求使用 delay_value 参数,该参数将创建网络参数的不可微分副本,用作目标网络。

loss_fn = DQNLoss(policy, action_space=env.action_spec, delay_value=True)

由于我们使用的是双 DQN,因此我们需要更新目标参数。我们将使用 SoftUpdate 实例来完成这项工作。

updater = SoftUpdate(loss_fn, eps=0.95)

optim = torch.optim.Adam(policy.parameters(), lr=3e-4)

收集器和回放缓冲区

我们构建了最简单的的数据收集器。我们将尝试用一百万帧训练我们的算法,每次扩展缓冲区 50 帧。缓冲区将被设计为存储 20000 个轨迹,每个轨迹 50 步。在每个优化步骤(每数据收集 16 个)中,我们将从缓冲区收集 4 个项目,总共 200 个转换。我们将使用 LazyMemmapStorage 存储将数据保存在磁盘上。

注意

为了提高效率,我们在这里只运行几千次迭代。在实际设置中,帧总数应设置为 1M。

collector = SyncDataCollector(env, stoch_policy, frames_per_batch=50, total_frames=200)
rb = TensorDictReplayBuffer(
    storage=LazyMemmapStorage(20_000), batch_size=4, prefetch=10
)

训练循环

为了跟踪进度,我们将在每次数据收集 50 次后在环境中运行一次策略,并在训练后绘制结果。

utd = 16
pbar = tqdm.tqdm(total=collector.total_frames)
longest = 0

traj_lens = []
for i, data in enumerate(collector):
    if i == 0:
        print(
            "Let us print the first batch of data.\nPay attention to the key names "
            "which will reflect what can be found in this data structure, in particular: "
            "the output of the QValueModule (action_values, action and chosen_action_value),"
            "the 'is_init' key that will tell us if a step is initial or not, and the "
            "recurrent_state keys.\n",
            data,
        )
    pbar.update(data.numel())
    # it is important to pass data that is not flattened
    rb.extend(data.unsqueeze(0).to_tensordict().cpu())
    for _ in range(utd):
        s = rb.sample().to(device, non_blocking=True)
        loss_vals = loss_fn(s)
        loss_vals["loss"].backward()
        optim.step()
        optim.zero_grad()
    longest = max(longest, data["step_count"].max().item())
    pbar.set_description(
        f"steps: {longest}, loss_val: {loss_vals['loss'].item(): 4.4f}, action_spread: {data['action'].sum(0)}"
    )
    exploration_module.step(data.numel())
    updater.step()

    with set_exploration_type(ExplorationType.MODE), torch.no_grad():
        rollout = env.rollout(10000, stoch_policy)
        traj_lens.append(rollout.get(("next", "step_count")).max().item())
  0%|          | 0/200 [00:00<?, ?it/s]Let us print the first batch of data.
Pay attention to the key names which will reflect what can be found in this data structure, in particular: the output of the QValueModule (action_values, action and chosen_action_value),the 'is_init' key that will tell us if a step is initial or not, and the recurrent_state keys.
 TensorDict(
    fields={
        action: Tensor(shape=torch.Size([50, 2]), device=cpu, dtype=torch.int64, is_shared=False),
        action_value: Tensor(shape=torch.Size([50, 2]), device=cpu, dtype=torch.float32, is_shared=False),
        chosen_action_value: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.float32, is_shared=False),
        collector: TensorDict(
            fields={
                traj_ids: Tensor(shape=torch.Size([50]), device=cpu, dtype=torch.int64, is_shared=False)},
            batch_size=torch.Size([50]),
            device=None,
            is_shared=False),
        done: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False),
        embed: Tensor(shape=torch.Size([50, 128]), device=cpu, dtype=torch.float32, is_shared=False),
        is_init: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False),
        next: TensorDict(
            fields={
                done: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False),
                is_init: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False),
                pixels: Tensor(shape=torch.Size([50, 1, 84, 84]), device=cpu, dtype=torch.float32, is_shared=False),
                recurrent_state_c: Tensor(shape=torch.Size([50, 1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
                recurrent_state_h: Tensor(shape=torch.Size([50, 1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
                reward: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.float32, is_shared=False),
                step_count: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.int64, is_shared=False),
                terminated: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False),
                truncated: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False)},
            batch_size=torch.Size([50]),
            device=None,
            is_shared=False),
        pixels: Tensor(shape=torch.Size([50, 1, 84, 84]), device=cpu, dtype=torch.float32, is_shared=False),
        recurrent_state_c: Tensor(shape=torch.Size([50, 1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
        recurrent_state_h: Tensor(shape=torch.Size([50, 1, 128]), device=cpu, dtype=torch.float32, is_shared=False),
        step_count: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.int64, is_shared=False),
        terminated: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False),
        truncated: Tensor(shape=torch.Size([50, 1]), device=cpu, dtype=torch.bool, is_shared=False)},
    batch_size=torch.Size([50]),
    device=None,
    is_shared=False)

 25%|██▌       | 50/200 [00:00<00:01, 124.16it/s]
 25%|██▌       | 50/200 [00:10<00:01, 124.16it/s]
steps: 11, loss_val:  0.0006, action_spread: tensor([ 4, 46]):  25%|██▌       | 50/200 [00:20<00:01, 124.16it/s]
steps: 11, loss_val:  0.0006, action_spread: tensor([ 4, 46]):  50%|█████     | 100/200 [00:20<00:24,  4.10it/s]
steps: 11, loss_val:  0.0006, action_spread: tensor([ 7, 43]):  50%|█████     | 100/200 [00:40<00:24,  4.10it/s]
steps: 11, loss_val:  0.0006, action_spread: tensor([ 7, 43]):  75%|███████▌  | 150/200 [00:41<00:15,  3.14it/s]
steps: 11, loss_val:  0.0006, action_spread: tensor([50,  0]):  75%|███████▌  | 150/200 [01:00<00:15,  3.14it/s]
steps: 11, loss_val:  0.0006, action_spread: tensor([50,  0]): 100%|██████████| 200/200 [01:01<00:00,  2.84it/s]
steps: 19, loss_val:  0.0005, action_spread: tensor([25, 25]): 100%|██████████| 200/200 [01:21<00:00,  2.84it/s]

让我们绘制我们的结果。

if traj_lens:
    from matplotlib import pyplot as plt

    plt.plot(traj_lens)
    plt.xlabel("Test collection")
    plt.title("Test trajectory lengths")
Test trajectory lengths

结论

我们已经了解了如何在 TorchRL 中将 RNN 集成到策略中。您现在应该能够

  • 创建一个充当 TensorDictModule 的 LSTM 模块。

  • 通过 InitTracker 变换指示 LSTM 模块需要重置。

  • 将此模块集成到策略和损失模块中。

  • 确保收集器知道循环状态条目,以便它们可以与其余数据一起存储在回放缓冲区中。

进一步阅读

  • TorchRL 文档可以在这里找到 here

脚本的总运行时间:(2 分 20.389 秒)

估计内存使用量:2212 MB

由 Sphinx-Gallery 生成的图库

文档

访问 PyTorch 的全面开发者文档

查看文档

教程

获取适合初学者和高级开发人员的深入教程

查看教程

资源

查找开发资源并获得问题的解答

查看资源